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

4
5
from typing import Any

6
from vllm.logger import init_logger
7
from vllm.model_executor.layers.batch_invariant import vllm_is_batch_invariant
8
from vllm.platforms import current_platform
9
10
11

logger = init_logger(__name__)

12
13
14
15
16
17
# Track whether upstream flash-attn is available on ROCm.
# Set during module initialization and never modified afterwards.
# This module-level flag avoids repeated import attempts and ensures
# consistent behavior (similar to IS_AITER_FOUND in _aiter_ops.py).
_ROCM_FLASH_ATTN_AVAILABLE = False

18
if current_platform.is_cuda():
19
    from vllm._custom_ops import reshape_and_cache_flash
20
21
22
23
24
    from vllm.vllm_flash_attn import (  # type: ignore[attr-defined]
        flash_attn_varlen_func,
        get_scheduler_metadata,
    )

25
elif current_platform.is_xpu():
26
    from vllm import _custom_ops as ops
27
    from vllm._xpu_ops import xpu_ops
28
29

    reshape_and_cache_flash = ops.reshape_and_cache_flash
30
31
    flash_attn_varlen_func = xpu_ops.flash_attn_varlen_func  # type: ignore[assignment]
    get_scheduler_metadata = xpu_ops.get_scheduler_metadata  # type: ignore[assignment]
32
33
elif current_platform.is_rocm():
    try:
34
        from flash_attn import flash_attn_varlen_func  # type: ignore[no-redef]
35
36
37

        # Mark that upstream flash-attn is available on ROCm
        _ROCM_FLASH_ATTN_AVAILABLE = True
38
39
    except ImportError:

40
        def flash_attn_varlen_func(*args: Any, **kwargs: Any) -> Any:  # type: ignore[no-redef,misc]
41
42
43
44
            raise ImportError(
                "ROCm platform requires upstream flash-attn "
                "to be installed. Please install flash-attn first."
            )
45

46
47
48
49
50
51
52
53
54
    # ROCm doesn't use scheduler metadata (FA3 feature), provide stub
    def get_scheduler_metadata(*args: Any, **kwargs: Any) -> None:  # type: ignore[misc]
        return None

    # ROCm uses the C++ custom op for reshape_and_cache
    from vllm import _custom_ops as ops

    reshape_and_cache_flash = ops.reshape_and_cache_flash

55

56
57
58
def get_flash_attn_version(
    requires_alibi: bool = False, head_size: int | None = None
) -> int | None:
59
60
    if current_platform.is_xpu():
        return 2
61
62
63
    if current_platform.is_rocm():
        # ROCm doesn't use vllm_flash_attn; return None to skip fa_version arg
        return None
64
65
    try:
        from vllm.vllm_flash_attn.flash_attn_interface import (
66
67
68
69
            fa_version_unsupported_reason,
            is_fa_version_supported,
        )

70
71
72
73
74
        device_capability = current_platform.get_device_capability()

        assert device_capability is not None

        # 1. default version depending on platform
75
76
77
78
79
80
81
82
83
        if device_capability.major == 9 and is_fa_version_supported(3):
            # Hopper (SM90): prefer FA3
            fa_version = 3
        elif device_capability.major == 10 and is_fa_version_supported(4):
            # Blackwell (SM100+, restrict to SM100 for now): prefer FA4
            fa_version = 4
        else:
            # Fallback to FA2
            fa_version = 2
84

85
        # 2. override if passed by environment or config
86
        from vllm.config import get_current_vllm_config_or_none
87

88
89
90
91
92
        vllm_config = get_current_vllm_config_or_none()
        if (
            vllm_config is not None
            and vllm_config.attention_config.flash_attn_version is not None
        ):
93
            fa_version = vllm_config.attention_config.flash_attn_version
94
95

        # 3. fallback for unsupported combinations
96
        if device_capability.major >= 10 and fa_version == 3:
97
            logger.warning_once(
98
                "Cannot use FA version 3 on Blackwell platform, "
99
                "defaulting to FA version 4 if supported, otherwise FA2."
100
            )
101
            fa_version = 4 if is_fa_version_supported(4) else 2
102
103

        if requires_alibi and fa_version == 3:
104
105
106
            logger.warning_once(
                "Cannot use FA version 3 with ALiBi, defaulting to FA version 2."
            )
107
108
            fa_version = 2

109
110
111
112
113
114
        if requires_alibi and fa_version == 4:
            logger.warning_once(
                "Cannot use FA version 4 with ALiBi, defaulting to FA version 2."
            )
            fa_version = 2

115
116
117
118
119
120
121
122
123
124
        # FA4 currently uses batch-shape-dependent scheduling
        # heuristics on SM100+, which breaks batch invariance.
        if vllm_is_batch_invariant() and fa_version == 4:
            logger.warning_once(
                "Cannot use FA version 4 with batch invariance, "
                "defaulting to FA version 2.",
                scope="local",
            )
            fa_version = 2

125
126
127
        # FA4 on SM100 (Blackwell) has TMEM capacity limits that restrict
        # supported head dimensions.
        # See: https://github.com/Dao-AILab/flash-attention/issues/1959
128
129
        # Exception: hdim 192 is supported for MLA's diff-headdim case
        # (qk=192, v=128), added upstream in commits 1a15733e/1b36ab19.
130
131
132
133
134
        if (
            fa_version == 4
            and device_capability.major >= 10
            and head_size is not None
            and head_size > 128
135
            and head_size != 192
136
137
138
139
140
141
142
143
        ):
            logger.warning_once(
                "FA4 on Blackwell does not support head_size=%d due to TMEM "
                "capacity limits, defaulting to FA version 2.",
                head_size,
            )
            fa_version = 2

144
        if not is_fa_version_supported(fa_version):
145
146
147
148
149
            logger.error(
                "Cannot use FA version %d is not supported due to %s",
                fa_version,
                fa_version_unsupported_reason(fa_version),
            )
150
151
152
153
154

        assert is_fa_version_supported(fa_version)
        return fa_version
    except (ImportError, AssertionError):
        return None
155
156
157


def flash_attn_supports_fp8() -> bool:
158
159
    return (
        get_flash_attn_version() == 3
160
        and current_platform.is_device_capability_family(90)
161
    )
162
163


164
165
166
167
168
169
170
def flash_attn_supports_sinks() -> bool:
    if current_platform.is_xpu():
        return True
    else:
        return get_flash_attn_version() == 3


171
172
def flash_attn_supports_mla():
    from vllm.platforms import current_platform
173

174
175
176
    if current_platform.is_cuda():
        try:
            from vllm.vllm_flash_attn.flash_attn_interface import (
177
178
179
                is_fa_version_supported,
            )

180
181
182
            return is_fa_version_supported(
                3
            ) and current_platform.is_device_capability_family(90)
183
184
185
186

            # NOTE(Lucas): FA4 CuteDSL does NOT currently support MLA's non-standard
            # head dimensions (576 for qk, 512 for v) due to TMEM capacity limits.

187
188
189
190
191
        except (ImportError, AssertionError):
            pass
    return False


192
def is_flash_attn_varlen_func_available() -> bool:
193
194
195
196
197
198
199
    """Check if flash_attn_varlen_func is available.

    This function determines whether the flash_attn_varlen_func imported at module
    level is a working implementation or a stub.

    Platform-specific sources:
    - CUDA: vllm.vllm_flash_attn.flash_attn_varlen_func
200
    - XPU: xpu_ops.flash_attn_varlen_func
201
202
203
204
205
206
207
208
209
210
211
212
213
214
215
216
217
218
219
    - ROCm: upstream flash_attn.flash_attn_varlen_func (if available)

    Note: This is separate from the AITER flash attention backend (rocm_aiter_fa.py)
    which uses rocm_aiter_ops.flash_attn_varlen_func. The condition to use AITER is
    handled separately via _aiter_ops.is_aiter_found_and_supported().

    Returns:
        bool: True if a working flash_attn_varlen_func implementation is available.
    """
    if current_platform.is_cuda() or current_platform.is_xpu():
        # CUDA and XPU always have flash_attn_varlen_func available
        return True

    if current_platform.is_rocm():
        # Use the flag set during module import to check if
        # upstream flash-attn was successfully imported
        return _ROCM_FLASH_ATTN_AVAILABLE

    return False