layer.py 15.6 KB
Newer Older
1
# SPDX-License-Identifier: Apache-2.0
2
"""Attention layer."""
3
from typing import Any, Dict, List, Optional
4
5
6

import torch
import torch.nn as nn
7
import torch.nn.functional as F
8

9
import vllm.envs as envs
10
from vllm.attention import AttentionType
11
from vllm.attention.selector import backend_name_to_enum, get_attn_backend
12
from vllm.config import CacheConfig, get_current_vllm_config
13
from vllm.forward_context import ForwardContext, get_forward_context
14
from vllm.model_executor.layers.linear import UnquantizedLinearMethod
15
16
from vllm.model_executor.layers.quantization.base_config import (
    QuantizationConfig)
17
from vllm.model_executor.layers.quantization.kv_cache import BaseKVCacheMethod
18
from vllm.platforms import _Backend, current_platform
19
from vllm.utils import direct_register_custom_op
20
21
22
23
24
25
26
27
28
29
30
31
32
33
34
35
36
37
38
39
40


class Attention(nn.Module):
    """Attention layer.

    This class takes query, key, and value tensors as input. The input tensors
    can either contain prompt tokens or generation tokens.
    The class does the following:

    1. Store the input key and value tensors in the KV cache.
    2. Perform (multi-head/multi-query/grouped-query) attention.
    3. Return the output tensor.
    """

    def __init__(
        self,
        num_heads: int,
        head_size: int,
        scale: float,
        num_kv_heads: Optional[int] = None,
        alibi_slopes: Optional[List[float]] = None,
41
        cache_config: Optional[CacheConfig] = None,
42
        quant_config: Optional[QuantizationConfig] = None,
43
        blocksparse_params: Optional[Dict[str, Any]] = None,
44
        logits_soft_cap: Optional[float] = None,
45
        per_layer_sliding_window: Optional[int] = None,
46
        use_mla: bool = False,
47
        prefix: str = "",
48
        attn_type: str = AttentionType.DECODER,
49
        **extra_impl_args,
50
    ) -> None:
51
52
53
54
        """
        The KV cache is stored inside this class and is accessed via
        `self.kv_cache`.
        """
55
        super().__init__()
56
57
58
59
60
61
62
63
64
        if per_layer_sliding_window is not None:
            # per-layer sliding window
            sliding_window = per_layer_sliding_window
        elif cache_config is not None:
            # model-level sliding window
            sliding_window = cache_config.sliding_window
        else:
            sliding_window = None

65
66
67
        if cache_config is not None:
            kv_cache_dtype = cache_config.cache_dtype
            block_size = cache_config.block_size
68
            is_attention_free = cache_config.is_attention_free
69
            calculate_kv_scales = cache_config.calculate_kv_scales
70
71
72
        else:
            kv_cache_dtype = "auto"
            block_size = 16
73
            is_attention_free = False
74
            calculate_kv_scales = False
75
76
        if num_kv_heads is None:
            num_kv_heads = num_heads
77

78
        # The default k/v_scale is set to 1.0. This is ignored
79
80
        # when kv-cache is not fp8, and should be used with
        # kv-cache in fp8_e5m2. For kv-cache in fp8_e4m3, we
81
        # expect the pre-quantized k/v_scale to be loaded along
82
83
        # with the model weights.
        self.kv_cache_dtype = kv_cache_dtype
84
85
86
87
88
89
90
91
92
        self.calculate_kv_scales = calculate_kv_scales
        self._k_scale = torch.tensor(1.0, dtype=torch.float32)
        self._v_scale = torch.tensor(1.0, dtype=torch.float32)

        # We also keep the float32 versions of k/v_scale for attention
        # backends that don't support tensors (Flashinfer)
        self._k_scale_float = 1.0
        self._v_scale_float = 1.0

93
        self.use_mla = use_mla
94
95
96
97
98
        self.num_heads = num_heads
        self.head_size = head_size
        self.num_kv_heads = num_kv_heads
        self.sliding_window = sliding_window

99
        quant_method = quant_config.get_quant_method(
100
            self, prefix=prefix) if quant_config else None
101
102
        if quant_method is not None and not isinstance(
                quant_method, UnquantizedLinearMethod):
103
            assert isinstance(quant_method, BaseKVCacheMethod)
104
105
            # TODO (mgoin): kv cache dtype should be specified in the FP8
            # checkpoint config and become the "auto" behavior
106
107
108
109
110
111
112
113
114
            if self.kv_cache_dtype == "fp8_e5m2":
                raise ValueError("fp8_e5m2 kv-cache is not supported with "
                                 "fp8 checkpoints.")
            # If quantization is enabled, we make "k_scale" and "v_scale"
            # parameters so that it can be loaded from the model checkpoint.
            # The k/v_scale will then be converted back to native float32
            # values after weight loading.
            self.quant_method = quant_method
            self.quant_method.create_weights(self)
115

116
117
118
        # During model initialization, the default dtype is set as the model
        # weight and activation dtype.
        dtype = torch.get_default_dtype()
119
120
121
122
123
124
125
        attn_backend = get_attn_backend(head_size,
                                        dtype,
                                        kv_cache_dtype,
                                        block_size,
                                        is_attention_free,
                                        blocksparse_params is not None,
                                        use_mla=use_mla)
126
        impl_cls = attn_backend.get_impl_cls()
127
        self.impl = impl_cls(num_heads, head_size, scale, num_kv_heads,
128
                             alibi_slopes, sliding_window, kv_cache_dtype,
129
130
                             blocksparse_params, logits_soft_cap, attn_type,
                             **extra_impl_args)
131
        self.backend = backend_name_to_enum(attn_backend.get_name())
132
        self.dtype = dtype
133

134
135
136
137
        # For cuda-alike (CUDA and ROCM) and cpu platforms, we control how
        # torch.compile works by registering the attention as one giant
        # opaque custom op. For other platforms, we directly call them
        # and let torch.compile handle them.
138
139
140
        self.use_direct_call = not current_platform.is_cuda_alike(
        ) and not current_platform.is_cpu()

141
        self.use_output = attn_backend.accept_output_buffer
142
143
144
145
146
        compilation_config = get_current_vllm_config().compilation_config
        if prefix in compilation_config.static_forward_context:
            raise ValueError(f"Duplicate layer name: {prefix}")
        compilation_config.static_forward_context[prefix] = self
        self.layer_name = prefix
147
        self.attn_type = attn_type
148
149
150
151
152
153
154
        # use a placeholder kv cache tensor during init, which will be replaced
        # by bind_kv_cache
        # this variable will not be accessed if use_direct_call is True
        self.kv_cache = [
            torch.tensor([]) for _ in range(get_current_vllm_config(
            ).parallel_config.pipeline_parallel_size)
        ]
155

156
157
158
        self.k_range = torch.tensor(envs.K_SCALE_CONSTANT, dtype=torch.float32)
        self.v_range = torch.tensor(envs.V_SCALE_CONSTANT, dtype=torch.float32)

159
160
161
162
163
    def forward(
        self,
        query: torch.Tensor,
        key: torch.Tensor,
        value: torch.Tensor,
164
165
166
167
        # For some alternate attention backends like MLA the attention output
        # shape does not match the query shape, so we optionally let the model
        # definition specify the output tensor shape.
        output_shape: Optional[torch.Size] = None,
168
    ) -> torch.Tensor:
169
170
171
172
173
174
175
176
177
        """
        The KV cache is stored inside this class and is accessed via
        `self.kv_cache`.

        Attention metadata (`attn_metadata`) is set using a context manager in
        the model runner's `execute_model` method. It is accessed via forward
        context using
        `vllm.forward_context.get_forward_context().attn_metadata`.
        """
Chen Zhang's avatar
Chen Zhang committed
178
        if self.calculate_kv_scales:
179
180
            attn_metadata = get_forward_context().attn_metadata
            if attn_metadata.enable_kv_scales_calculation:
Chen Zhang's avatar
Chen Zhang committed
181
                self.calc_kv_scales(key, value)
182
        if self.use_output:
183
184
185
186
187
188
189
190
191
192
193
194
195
196
197
198
199
200
201
            output_shape = (output_shape
                            if output_shape is not None else query.shape)
            output = torch.empty(output_shape,
                                 dtype=query.dtype,
                                 device=query.device)
            hidden_size = output_shape[-1]
            # We skip reshaping query, key and value tensors for the MLA
            # backend since these tensors have different semantics and are
            # processed differently.
            if not self.use_mla:
                # Reshape the query, key, and value tensors.
                # NOTE(woosuk): We do this outside the custom op to minimize the
                # CPU overheads from the non-CUDA-graph regions.
                query = query.view(-1, self.num_heads, self.head_size)
                output = output.view(-1, self.num_heads, self.head_size)
                if key is not None:
                    key = key.view(-1, self.num_kv_heads, self.head_size)
                if value is not None:
                    value = value.view(-1, self.num_kv_heads, self.head_size)
202
            if self.use_direct_call:
Chen Zhang's avatar
Chen Zhang committed
203
                forward_context: ForwardContext = get_forward_context()
204
                attn_metadata = forward_context.attn_metadata
Chen Zhang's avatar
Chen Zhang committed
205
206
207
208
209
210
                self_kv_cache = self.kv_cache[forward_context.virtual_engine]
                self.impl.forward(self,
                                  query,
                                  key,
                                  value,
                                  self_kv_cache,
211
                                  attn_metadata,
Chen Zhang's avatar
Chen Zhang committed
212
                                  output=output)
213
214
215
            else:
                torch.ops.vllm.unified_attention_with_output(
                    query, key, value, output, self.layer_name)
216
            return output.view(-1, hidden_size)
217
        else:
218
            if self.use_direct_call:
Chen Zhang's avatar
Chen Zhang committed
219
                forward_context = get_forward_context()
220
                attn_metadata = forward_context.attn_metadata
Chen Zhang's avatar
Chen Zhang committed
221
222
                self_kv_cache = self.kv_cache[forward_context.virtual_engine]
                return self.impl.forward(self, query, key, value,
223
                                         self_kv_cache, attn_metadata)
224
225
226
            else:
                return torch.ops.vllm.unified_attention(
                    query, key, value, self.layer_name)
227

228
229
230
231
232
233
234
235
    def calc_kv_scales(self, key, value):
        self._k_scale.copy_(torch.abs(key).max() / self.k_range)
        self._v_scale.copy_(torch.abs(value).max() / self.v_range)
        self._k_scale_float = self._k_scale.item()
        self._v_scale_float = self._v_scale.item()
        # We only calculate the scales once
        self.calculate_kv_scales = False

236
237
238
239
240
    def extra_repr(self) -> str:
        s = f"head_size={self.impl.head_size}"  # type: ignore
        s += f", num_heads={self.impl.num_heads}"  # type: ignore
        s += f", num_kv_heads={self.impl.num_kv_heads}"  # type: ignore
        s += f", scale={self.impl.scale}"  # type: ignore
241
        s += f", backend={self.impl.__class__.__name__}"
242
        return s
243

244
    def process_weights_after_loading(self, act_dtype: torch.dtype):
245
        if hasattr(self.impl, "process_weights_after_loading"):
246
            self.impl.process_weights_after_loading(act_dtype)
247

248

249
250
251
252
253
254
255
256
257
258
259
260
261
262
263
264
class MultiHeadAttention(nn.Module):
    """Multi-headed attention without any cache, used for ViT."""

    def __init__(
        self,
        num_heads: int,
        head_size: int,
        scale: float,
        num_kv_heads: Optional[int] = None,
    ):
        super().__init__()
        self.num_heads = num_heads
        self.head_size = head_size
        self.scale = scale
        self.num_kv_heads = num_heads if num_kv_heads is None else num_kv_heads

265
266
267
        assert self.num_heads % self.num_kv_heads == 0
        self.num_queries_per_kv = self.num_heads // self.num_kv_heads

268
269
270
271
272
273
        dtype = torch.get_default_dtype()
        attn_backend = get_attn_backend(head_size,
                                        dtype,
                                        kv_cache_dtype=None,
                                        block_size=16,
                                        is_attention_free=False)
274
        backend = backend_name_to_enum(attn_backend.get_name())
275
276
        if backend in {_Backend.FLASH_ATTN, _Backend.FLASH_ATTN_VLLM_V1}:
            backend = _Backend.XFORMERS
277

278
        self.attn_backend = backend if backend in {
279
280
            _Backend.TORCH_SDPA,
            _Backend.XFORMERS,
281
282
283
284
285
286
287
288
289
        } else _Backend.TORCH_SDPA

    def forward(
        self,
        query: torch.Tensor,
        key: torch.Tensor,
        value: torch.Tensor,
    ) -> torch.Tensor:
        """Input shape: batch_size x seq_len x hidden_size"""
290
        # TODO(Isotr0py): Use existing backend implementations and support FA3
291
292
293
294
295
296
297
        bsz, q_len, _ = query.size()
        kv_len = key.size(1)

        query = query.view(bsz, q_len, self.num_heads, self.head_size)
        key = key.view(bsz, kv_len, self.num_kv_heads, self.head_size)
        value = value.view(bsz, kv_len, self.num_kv_heads, self.head_size)

298
299
300
301
302
        if (num_repeat := self.num_queries_per_kv) > 1:
            # Handle MQA and GQA
            key = torch.repeat_interleave(key, num_repeat, dim=2)
            value = torch.repeat_interleave(value, num_repeat, dim=2)

303
        if self.attn_backend == _Backend.XFORMERS:
304
305
306
307
308
309
310
311
312
313
314
315
316
317
            from xformers import ops as xops

            out = xops.memory_efficient_attention_forward(query,
                                                          key,
                                                          value,
                                                          scale=self.scale)
        elif self.attn_backend == _Backend.TORCH_SDPA:
            query, key, value = (x.transpose(1, 2)
                                 for x in (query, key, value))
            out = F.scaled_dot_product_attention(query,
                                                 key,
                                                 value,
                                                 scale=self.scale)
            out = out.transpose(1, 2)
318
        return out.reshape(bsz, q_len, -1)
319
320


321
322
323
324
325
326
327
def unified_attention(
    query: torch.Tensor,
    key: torch.Tensor,
    value: torch.Tensor,
    layer_name: str,
) -> torch.Tensor:
    forward_context: ForwardContext = get_forward_context()
328
    attn_metadata = forward_context.attn_metadata
329
    self = forward_context.no_compile_layers[layer_name]
330
    kv_cache = self.kv_cache[forward_context.virtual_engine]
331
    return self.impl.forward(self, query, key, value, kv_cache, attn_metadata)
332
333
334
335
336
337
338
339
340
341
342
343
344
345


def unified_attention_fake(
    query: torch.Tensor,
    key: torch.Tensor,
    value: torch.Tensor,
    layer_name: str,
) -> torch.Tensor:
    return torch.empty_like(query).contiguous()


direct_register_custom_op(
    op_name="unified_attention",
    op_func=unified_attention,
346
    mutates_args=[],
347
348
349
    fake_impl=unified_attention_fake,
    dispatch_key=current_platform.dispatch_key,
)
350
351
352
353
354
355
356
357
358
359


def unified_attention_with_output(
    query: torch.Tensor,
    key: torch.Tensor,
    value: torch.Tensor,
    output: torch.Tensor,
    layer_name: str,
) -> None:
    forward_context: ForwardContext = get_forward_context()
360
    attn_metadata = forward_context.attn_metadata
361
    self = forward_context.no_compile_layers[layer_name]
362
    kv_cache = self.kv_cache[forward_context.virtual_engine]
363
364
    self.impl.forward(self,
                      query,
365
366
367
368
369
370
371
372
373
374
375
376
377
378
379
380
381
382
383
384
                      key,
                      value,
                      kv_cache,
                      attn_metadata,
                      output=output)


def unified_attention_with_output_fake(
    query: torch.Tensor,
    key: torch.Tensor,
    value: torch.Tensor,
    output: torch.Tensor,
    layer_name: str,
) -> None:
    return


direct_register_custom_op(
    op_name="unified_attention_with_output",
    op_func=unified_attention_with_output,
385
    mutates_args=["output"],
386
387
388
    fake_impl=unified_attention_with_output_fake,
    dispatch_key=current_platform.dispatch_key,
)