flashmla.py 6.54 KB
Newer Older
1
# SPDX-License-Identifier: Apache-2.0
2
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
3
4

from dataclasses import dataclass
5
from typing import ClassVar, Optional
6
7
8

import torch

9
10
from vllm.attention.backends.abstract import (AttentionType,
                                              is_quantized_kv_cache)
11
12
13
from vllm.attention.ops.flashmla import (flash_mla_with_kvcache,
                                         get_mla_metadata,
                                         is_flashmla_supported)
14
from vllm.config import VllmConfig
15
16
from vllm.logger import init_logger
from vllm.v1.attention.backends.mla.common import (MLACommonBackend,
17
                                                   MLACommonDecodeMetadata,
18
19
20
                                                   MLACommonImpl,
                                                   MLACommonMetadata,
                                                   MLACommonMetadataBuilder)
21
from vllm.v1.kv_cache_interface import AttentionSpec
22
23
24
25
26
27
28
29
30
31
32

logger = init_logger(__name__)


class FlashMLABackend(MLACommonBackend):

    @staticmethod
    def get_name() -> str:
        return "FLASHMLA_VLLM_V1"

    @staticmethod
33
    def get_metadata_cls() -> type["FlashMLAMetadata"]:
34
35
36
        return FlashMLAMetadata

    @staticmethod
37
    def get_builder_cls() -> type["FlashMLAMetadataBuilder"]:
38
39
40
        return FlashMLAMetadataBuilder

    @staticmethod
41
    def get_impl_cls() -> type["FlashMLAImpl"]:
42
43
44
45
        return FlashMLAImpl


@dataclass
46
class FlashMLADecodeMetadata(MLACommonDecodeMetadata):
47
    tile_scheduler_metadata: torch.Tensor
48
49
50
51
52
53
    num_splits: torch.Tensor


@dataclass
class FlashMLAMetadata(MLACommonMetadata[FlashMLADecodeMetadata]):
    pass
54
55
56


class FlashMLAMetadataBuilder(MLACommonMetadataBuilder[FlashMLAMetadata]):
57
    full_cudagraph_supported: ClassVar[bool] = True  # Decode-only
58

59
60
61
    def __init__(self, kv_cache_spec: AttentionSpec, vllm_config: VllmConfig,
                 device: torch.device):
        super().__init__(kv_cache_spec, vllm_config, device, FlashMLAMetadata)
62

63
64
65
        self.compilation_config = vllm_config.compilation_config
        self.num_q_heads = vllm_config.model_config.get_num_attention_heads(
            vllm_config.parallel_config)
66

67
68
69
        self.cg_buf_tile_scheduler_metadata = None
        self.cg_buf_num_splits = None

70
    def _build_decode(self, block_table_tensor: torch.Tensor,
71
72
73
74
75
76
77
                      seq_lens: torch.Tensor) -> FlashMLADecodeMetadata:
        tile_scheduler_metadata, num_splits = \
            get_mla_metadata(
            seq_lens,
            self.num_q_heads,
            1, # MQA for the decode path
        )
78

79
        if self.compilation_config.full_cuda_graph:
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
            # First time around (CUDAGraph capture), allocate the static buffer
            if self.cg_buf_tile_scheduler_metadata is None:
                self.cg_buf_tile_scheduler_metadata = tile_scheduler_metadata
                self.cg_buf_num_splits = num_splits
            else:
                assert self.cg_buf_num_splits is not None

                # Metadata per-SM, fixed size (#SMs, TileMetadataSize)
                assert (self.cg_buf_tile_scheduler_metadata.size() ==
                        tile_scheduler_metadata.size())
                self.cg_buf_tile_scheduler_metadata.\
                    copy_(tile_scheduler_metadata)
                tile_scheduler_metadata = self.cg_buf_tile_scheduler_metadata

                # Num splits is per-batch, varying size (batch_size,)
                n = num_splits.size(0)
                # make sure static buffer is large enough
                assert n <= self.cg_buf_num_splits.size(0)
                num_splits_view = self.cg_buf_num_splits[:n]
                num_splits_view.copy_(num_splits)
                self.cg_buf_num_splits[n:].fill_(0)  # fill the rest with 0s
                num_splits = num_splits_view

103
        return FlashMLADecodeMetadata(
104
            block_table=block_table_tensor,
105
106
107
108
            seq_lens=seq_lens,
            tile_scheduler_metadata=tile_scheduler_metadata,
            num_splits=num_splits,
        )
109
110
111
112
113
114
115
116
117
118


class FlashMLAImpl(MLACommonImpl[FlashMLAMetadata]):

    def __init__(
            self,
            num_heads: int,
            head_size: int,
            scale: float,
            num_kv_heads: int,
119
            alibi_slopes: Optional[list[float]],
120
121
122
123
            sliding_window: Optional[int],
            kv_cache_dtype: str,
            logits_soft_cap: Optional[float],
            attn_type: str,
124
            kv_sharing_target_layer_name: Optional[str],
125
126
127
128
            # MLA Specific Arguments
            **mla_args) -> None:
        super().__init__(num_heads, head_size, scale, num_kv_heads,
                         alibi_slopes, sliding_window, kv_cache_dtype,
129
                         logits_soft_cap, attn_type,
130
                         kv_sharing_target_layer_name, **mla_args)
131
132
133
134

        assert is_flashmla_supported(), \
            "FlashMLA is not supported on this device"

135
        unsupported_features = [alibi_slopes, sliding_window, logits_soft_cap]
136
137
138
        if any(unsupported_features):
            raise NotImplementedError(
                "FlashMLAImpl does not support one of the following: "
139
                "alibi_slopes, sliding_window, logits_soft_cap")
140
141
142
143
144
145
146

        if attn_type != AttentionType.DECODER:
            raise NotImplementedError("Encoder self-attention and "
                                      "encoder/decoder cross-attention "
                                      "are not implemented for "
                                      "FlashMLAImpl")

147
148
149
150
        if is_quantized_kv_cache(self.kv_cache_dtype):
            raise NotImplementedError(
                "FlashMLA V1 with FP8 KV cache not yet supported")

151
152
153
154
155
156
157
158
    def _forward_decode(
        self,
        q_nope: torch.Tensor,
        q_pe: torch.Tensor,
        kv_c_and_k_pe_cache: torch.Tensor,
        attn_metadata: FlashMLAMetadata,
    ) -> torch.Tensor:
        assert kv_c_and_k_pe_cache.numel() > 0
159
160
        assert attn_metadata.decode is not None

161
162
163
164
165
166
        q = torch.cat([q_nope, q_pe], dim=-1)\
            .unsqueeze(1) # Add seqlen dim of 1 (decode)

        o, _ = flash_mla_with_kvcache(
            q=q,
            k_cache=kv_c_and_k_pe_cache.unsqueeze(-2),  # Add head dim of 1
167
168
            block_table=attn_metadata.decode.block_table,
            cache_seqlens=attn_metadata.decode.seq_lens,
169
            head_dim_v=self.kv_lora_rank,
170
171
172
            tile_scheduler_metadata=attn_metadata.decode.
            tile_scheduler_metadata,
            num_splits=attn_metadata.decode.num_splits,
173
174
175
176
            softmax_scale=self.scale,
            causal=True,
        )

177
        return self._v_up_proj(o)