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

4
from unittest.mock import patch
5
6
7
8

import pytest
import torch

9
from vllm.attention.selector import _cached_get_attn_backend, get_attn_backend
10
from vllm.platforms import current_platform
11
12
13
from vllm.platforms.cpu import CpuPlatform
from vllm.platforms.cuda import CudaPlatform
from vllm.platforms.rocm import RocmPlatform
14
from vllm.utils import STR_BACKEND_ENV_VAR, STR_FLASH_ATTN_VAL, STR_INVALID_VAL
15
16


17
18
@pytest.fixture(autouse=True)
def clear_cache():
19
    """Clear lru cache to ensure each test case runs without caching."""
20
21
22
    _cached_get_attn_backend.cache_clear()


23
24
# Define MLA and non-MLA backends separately
DEVICE_MLA_BACKENDS = {
25
    "cuda": [
26
27
28
29
30
        "TRITON_MLA",
        "FLASHMLA",
        "FLASHINFER_MLA",
        "FLASH_ATTN_MLA",
        "CUTLASS_MLA",
31
    ],
32
33
34
35
36
    "hip": ["TRITON_MLA", "ROCM_AITER_MLA"],
    "cpu": [],
}

DEVICE_REGULAR_ATTN_BACKENDS = {
37
    "cuda": ["XFORMERS", "FLASHINFER", "FLASH_ATTN"],
38
    "hip": ["ROCM_ATTN"],
39
    "cpu": ["CPU_ATTN"],
40
41
42
43
44
}

DEVICE_MLA_BLOCK_SIZES = {
    "cuda": [16, 64],  # CUDA supports both standard and extended block sizes
    "hip": [16, 1],  # HIP requires special handling for block_size=1
45
    # "cpu": [16]  # CPU uses fixed block size from test cases
46
    "cpu": [],  # FIXME(woosuk): Temporarily disable CPU tests
47
48
49
50
}


def generate_params():
51
    is_rocm = current_platform.is_rocm()
52
    params = []
53
    device_list = ["cuda", "cpu"] if not is_rocm else ["hip", "cpu"]
54
    for use_mla in [True, False]:
55
        for device in device_list:
56
57
58
59
60
            backends = (
                DEVICE_MLA_BACKENDS[device]
                if use_mla
                else DEVICE_REGULAR_ATTN_BACKENDS[device]
            )
61
            for name in backends:
62
                block_sizes = DEVICE_MLA_BLOCK_SIZES[device] if use_mla else [16]
63
64
65
66
67
68
69
                for block_size in block_sizes:
                    params.append(
                        pytest.param(
                            device,
                            name,
                            use_mla,
                            block_size,
70
71
72
                            id=f"{device}_{name}_mla_{str(use_mla)[0]}_blks{block_size}",
                        )
                    )
73
74
75
    return params


76
@pytest.mark.parametrize("device, name, use_mla, block_size", generate_params())
77
def test_env(
78
    device: str,
79
    name: str,
80
81
    use_mla: bool,
    block_size: int,
82
83
    monkeypatch: pytest.MonkeyPatch,
):
84
    """Test attention backend selection with valid device-backend pairs."""
85
86
    with monkeypatch.context() as m:
        m.setenv(STR_BACKEND_ENV_VAR, name)
87
        m.setenv("VLLM_MLA_DISABLE", "1" if use_mla else "0")
88
89

        if device == "cpu":
90
            with patch("vllm.platforms.current_platform", CpuPlatform()):
91
                backend = get_attn_backend(16, torch.float16, None, block_size)
92
            assert backend.get_name() == "CPU_ATTN"
93

94
        elif device == "hip":
95
            with patch("vllm.platforms.current_platform", RocmPlatform()):
96
                if use_mla:
97
98
99
100
101
102
103
104
                    # ROCm MLA backend logic:
                    # - TRITON_MLA: supported when block_size != 1
                    # - ROCM_AITER_MLA: supported when block_size == 1
                    # If backend is forced but doesn't match block_size,
                    # should raise ValueError

                    if name == "TRITON_MLA" and block_size == 1:
                        # TRITON_MLA doesn't support block_size == 1
105
                        with pytest.raises(ValueError) as exc_info:
106
107
108
109
                            get_attn_backend(
                                16, torch.float16, None, block_size, use_mla=use_mla
                            )
                        assert f"The selected backend, {name}" in str(exc_info.value)
110
111
                    else:
                        # Valid backend-block_size combination
112
113
114
                        backend = get_attn_backend(
                            16, torch.float16, None, block_size, use_mla=use_mla
                        )
115
                        expected = name
116
                        assert backend.get_name() == expected
117
                else:
118
119
120
                    backend = get_attn_backend(
                        16, torch.float16, None, block_size, use_mla=use_mla
                    )
121
                    expected = "ROCM_ATTN"
122
123
124
                    assert backend.get_name() == expected

        elif device == "cuda":
125
            with patch("vllm.platforms.current_platform", CudaPlatform()):
126
                capability = torch.cuda.get_device_capability()
127
                if use_mla:
128
129
                    # CUDA MLA backend logic:
                    # - CUTLASS_MLA: only supported with block_size == 128
130
                    #   and Blackwell GPUs (SM 10.x), V1 only
131
                    # - FLASHINFER_MLA: only supported on Blackwell GPUs
132
                    #   (SM 10.x), V1 only
133
134
135
136
137
                    # - FLASHMLA: only supported with block_size == 64
                    # - FLASH_ATTN_MLA: V1 only
                    # - TRITON_MLA: fallback for other cases

                    if name == "CUTLASS_MLA":
138
                        if block_size != 128:
139
                            # CUTLASS_MLA only supports block_size == 128
140
                            pytest.skip("CUTLASS_MLA only supports block_size 128")
141
142
143
144
145
146
147
                        if capability[0] != 10:
                            pytest.skip("CUTLASS MLA is not supported on this platform")
                        backend = get_attn_backend(
                            576, torch.float16, None, block_size, use_mla=use_mla
                        )
                        expected = "CUTLASS_MLA"
                        assert backend.get_name() == expected
148
                    elif name == "FLASHINFER_MLA":
149
150
151
152
                        if capability[0] != 10:
                            pytest.skip(
                                "FlashInfer MLA is not supported on this platform"
                            )
153
                        if block_size not in [32, 64]:
154
155
                            # FlashInfer MLA only supports block_size 32 or 64
                            pytest.skip(
156
157
                                "FlashInfer MLA only supports block_size 32 or 64"
                            )
158
159
160
161
162
                        backend = get_attn_backend(
                            576, torch.float16, None, block_size, use_mla=use_mla
                        )
                        expected = "FLASHINFER_MLA"
                        assert backend.get_name() == expected
163
164
165
166
                    elif name == "FLASHMLA":
                        if block_size != 64:
                            # FlashMLA only supports block_size == 64
                            pytest.skip("FlashMLA only supports block_size 64")
167
168
169
                        from vllm.v1.attention.backends.mla.flashmla import (
                            is_flashmla_dense_supported,
                        )
170

171
172
173
174
175
176
177
178
179
180
181
182
                        is_supported, _ = is_flashmla_dense_supported()
                        if not is_supported:
                            pytest.skip("FlashMLA not supported on this platform")
                        backend = get_attn_backend(
                            576,
                            torch.float16,
                            None,
                            block_size,
                            use_mla=use_mla,
                        )
                        expected = name
                        assert backend.get_name() == expected
183
                    elif name == "FLASH_ATTN_MLA":
184
185
186
187
188
189
190
191
                        from vllm.attention.utils.fa_utils import (
                            flash_attn_supports_mla,
                        )

                        if not flash_attn_supports_mla():
                            pytest.skip(
                                "FlashAttention MLA not supported on this platform"
                            )
192
                        backend = get_attn_backend(
193
                            576, torch.float16, None, block_size, use_mla=use_mla
194
                        )
195
196
                        expected = "FLASH_ATTN_MLA"
                        assert backend.get_name() == expected
197
                    else:
198
                        # TRITON_MLA or other fallback
199
                        backend = get_attn_backend(
200
                            576, torch.float16, None, block_size, use_mla=use_mla
201
                        )
202
                        expected = "TRITON_MLA"
203
                        assert backend.get_name() == expected
204
                elif name == "FLASHINFER":
205
                    backend = get_attn_backend(
206
                        64, torch.float16, None, block_size, use_mla=use_mla
207
                    )
208
                    expected = "FLASHINFER"
209
                    assert backend.get_name() == expected
210
                elif name == "XFORMERS":
211
212
213
                    backend = get_attn_backend(
                        32, torch.float16, None, block_size, use_mla=use_mla
                    )
214
                    expected = "XFORMERS"
215
                    assert backend.get_name() == expected
216
                elif name == "FLASH_ATTN":
217
218
219
                    backend = get_attn_backend(
                        32, torch.float16, None, block_size, use_mla=use_mla
                    )
220
221
                    expected = "FLASH_ATTN"
                    assert backend.get_name() == expected
222

223

224
@pytest.mark.parametrize("device", ["cpu", "cuda"])
225
def test_fp32_fallback(device: str):
226
    """Test attention backend selection with fp32."""
227
    if device == "cpu":
228
        with patch("vllm.platforms.current_platform", CpuPlatform()):
229
            backend = get_attn_backend(16, torch.float32, None, 16)
230
        assert backend.get_name() == "CPU_ATTN"
231

232
    elif device == "cuda":
233
        with patch("vllm.platforms.current_platform", CudaPlatform()):
234
235
            backend = get_attn_backend(16, torch.float32, None, 16)
        assert backend.get_name() == "FLEX_ATTENTION"
236
237


238
def test_flash_attn(monkeypatch: pytest.MonkeyPatch):
239
    """Test FlashAttn validation."""
240
241
242
243
    pytest.skip(
        "Skipping as current backend selector does not "
        "handle fallbacks when a backend is set via env var."
    )
244

245
246
    with monkeypatch.context() as m:
        m.setenv(STR_BACKEND_ENV_VAR, STR_FLASH_ATTN_VAL)
247

248
        # Unsupported CUDA arch
249
        monkeypatch.setattr(torch.cuda, "get_device_capability", lambda _=None: (7, 5))
250
        backend = get_attn_backend(16, torch.float16, None, 16)
251
        assert backend.get_name() != STR_FLASH_ATTN_VAL
252

253
254
        # Reset the monkeypatch for subsequent tests
        monkeypatch.undo()
255

256
        # Unsupported data type
257
        backend = get_attn_backend(16, torch.float8_e4m3fn, None, 16)
258
        assert backend.get_name() != STR_FLASH_ATTN_VAL
259

260
        # Unsupported kv cache data type
261
        backend = get_attn_backend(16, torch.float16, "fp8", 16)
262
        assert backend.get_name() != STR_FLASH_ATTN_VAL
263

264
        # Unsupported block size
265
        backend = get_attn_backend(16, torch.float16, None, 8)
266
267
268
269
        assert backend.get_name() != STR_FLASH_ATTN_VAL

        # flash-attn is not installed
        import sys
270
271
272

        original_module = sys.modules.get("vllm_flash_attn")
        monkeypatch.setitem(sys.modules, "vllm_flash_attn", None)
273
        backend = get_attn_backend(16, torch.float16, None, 16)
274
        assert backend.get_name() != STR_FLASH_ATTN_VAL
275

276
277
        # Restore the original module if it existed
        if original_module is not None:
278
            monkeypatch.setitem(sys.modules, "vllm_flash_attn", original_module)
279
        else:
280
            monkeypatch.delitem(sys.modules, "vllm_flash_attn", raising=False)
281

282
        # Unsupported head size
283
        backend = get_attn_backend(17, torch.float16, None, 16)
284
        assert backend.get_name() != STR_FLASH_ATTN_VAL
285
286


287
def test_invalid_env(monkeypatch: pytest.MonkeyPatch):
288
    """Test that invalid attention backend names raise ValueError."""
289
290
    with (
        monkeypatch.context() as m,
291
        patch("vllm.platforms.current_platform", CudaPlatform()),
292
    ):
293
        m.setenv(STR_BACKEND_ENV_VAR, STR_INVALID_VAL)
294

295
296
        # Should raise ValueError for invalid backend
        with pytest.raises(ValueError) as exc_info:
297
            get_attn_backend(32, torch.float16, None, 16)
298
        assert "Invalid value 'INVALID'" in str(exc_info.value)