fa_utils.py 4.1 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.platforms import current_platform
8
9
10

logger = init_logger(__name__)

11
if current_platform.is_cuda():
12
    from vllm._custom_ops import reshape_and_cache_flash
13
14
15
16
17
    from vllm.vllm_flash_attn import (  # type: ignore[attr-defined]
        flash_attn_varlen_func,
        get_scheduler_metadata,
    )

18
elif current_platform.is_xpu():
19
    from vllm._ipex_ops import ipex_ops
20

21
    reshape_and_cache_flash = ipex_ops.reshape_and_cache_flash
22
23
    flash_attn_varlen_func = ipex_ops.flash_attn_varlen_func  # type: ignore[assignment]
    get_scheduler_metadata = ipex_ops.get_scheduler_metadata  # type: ignore[assignment]
24

25
26
elif current_platform.is_rocm():
    try:
27
        from flash_attn import flash_attn_varlen_func  # type: ignore[no-redef]
28
29
    except ImportError:

30
        def flash_attn_varlen_func(*args: Any, **kwargs: Any) -> Any:  # type: ignore[no-redef,misc]
31
32
33
34
            raise ImportError(
                "ROCm platform requires upstream flash-attn "
                "to be installed. Please install flash-attn first."
            )
35

36

37
def get_flash_attn_version(requires_alibi: bool = False) -> int | None:
38
39
    # import here to avoid circular dependencies
    from vllm.platforms import current_platform
40

41
42
    if current_platform.is_xpu():
        return 2
43
44
45
    if current_platform.is_rocm():
        # ROCm doesn't use vllm_flash_attn; return None to skip fa_version arg
        return None
46
47
    try:
        from vllm.vllm_flash_attn.flash_attn_interface import (
48
49
50
51
            fa_version_unsupported_reason,
            is_fa_version_supported,
        )

52
53
54
55
56
        device_capability = current_platform.get_device_capability()

        assert device_capability is not None

        # 1. default version depending on platform
57
58
59
        fa_version = (
            3 if (device_capability.major == 9 and is_fa_version_supported(3)) else 2
        )
60

61
        # 2. override if passed by environment or config
62
        from vllm.config import get_current_vllm_config_or_none
63

64
65
66
67
68
        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
        ):
69
            fa_version = vllm_config.attention_config.flash_attn_version
70
71
72

        # 3. fallback for unsupported combinations
        if device_capability.major == 10 and fa_version == 3:
73
            logger.warning_once(
74
                "Cannot use FA version 3 on Blackwell platform, "
75
76
                "defaulting to FA version 2."
            )
77
78
79
            fa_version = 2

        if requires_alibi and fa_version == 3:
80
81
82
            logger.warning_once(
                "Cannot use FA version 3 with ALiBi, defaulting to FA version 2."
            )
83
84
85
            fa_version = 2

        if not is_fa_version_supported(fa_version):
86
87
88
89
90
            logger.error(
                "Cannot use FA version %d is not supported due to %s",
                fa_version,
                fa_version_unsupported_reason(fa_version),
            )
91
92
93
94
95

        assert is_fa_version_supported(fa_version)
        return fa_version
    except (ImportError, AssertionError):
        return None
96
97
98


def flash_attn_supports_fp8() -> bool:
99
100
    return (
        get_flash_attn_version() == 3
101
        and current_platform.is_device_capability_family(90)
102
    )
103
104


105
106
107
108
109
110
111
def flash_attn_supports_sinks() -> bool:
    if current_platform.is_xpu():
        return True
    else:
        return get_flash_attn_version() == 3


112
113
def flash_attn_supports_mla():
    from vllm.platforms import current_platform
114

115
116
117
    if current_platform.is_cuda():
        try:
            from vllm.vllm_flash_attn.flash_attn_interface import (
118
119
120
                is_fa_version_supported,
            )

121
122
123
            return is_fa_version_supported(
                3
            ) and current_platform.is_device_capability_family(90)
124
125
126
127
128
        except (ImportError, AssertionError):
            pass
    return False


129
130
def is_flash_attn_varlen_func_available() -> bool:
    return current_platform.is_cuda() or current_platform.is_xpu()