"mre/fastpt-overhead/scripts/run_with_probe.sh" did not exist on "63d618ba7bcc39fd245837ecc1f6a7954e079409"
rocm.py 17.7 KB
Newer Older
1
# SPDX-License-Identifier: Apache-2.0
2
# SPDX-FileCopyrightText: Copyright contributors to the vLLM project
3

4
import os
5
from functools import cache, lru_cache, wraps
6
from typing import TYPE_CHECKING
7
8
9

import torch

10
import vllm.envs as envs
11
from vllm.logger import init_logger
12
from vllm.utils import cuda_device_count_stateless
13

14
from .interface import DeviceCapability, Platform, PlatformEnum
15

16
if TYPE_CHECKING:
17
    from vllm.attention.backends.registry import _Backend
18
    from vllm.config import ModelConfig, VllmConfig
19
20
else:
    _Backend = None
21

22
23
logger = init_logger(__name__)

24
try:
25
26
27
28
29
30
31
32
    from amdsmi import (
        AmdSmiException,
        amdsmi_get_gpu_asic_info,
        amdsmi_get_processor_handles,
        amdsmi_init,
        amdsmi_shut_down,
        amdsmi_topo_get_link_type,
    )
33
34
35
except ImportError as e:
    logger.warning("Failed to import from amdsmi with %r", e)

36
37
38
39
40
41
42
43
44
45
46
try:
    import vllm._C  # noqa: F401
except ImportError as e:
    logger.warning("Failed to import from vllm._C with %r", e)

# import custom ops, trigger op registration
try:
    import vllm._rocm_C  # noqa: F401
except ImportError as e:
    logger.warning("Failed to import from vllm._rocm_C with %r", e)

47
# Models not supported by ROCm.
48
_ROCM_UNSUPPORTED_MODELS: list[str] = []
49
50
51

# Models partially supported by ROCm.
# Architecture -> Reason.
52
53
54
55
56
57
_ROCM_SWA_REASON = (
    "Sliding window attention (SWA) is not yet supported in "
    "Triton flash attention. For half-precision SWA support, "
    "please use CK flash attention by setting "
    "`VLLM_USE_TRITON_FLASH_ATTN=0`"
)
58
_ROCM_PARTIALLY_SUPPORTED_MODELS: dict[str, str] = {
59
60
61
62
63
64
65
66
    "Qwen2ForCausalLM": _ROCM_SWA_REASON,
    "MistralForCausalLM": _ROCM_SWA_REASON,
    "MixtralForCausalLM": _ROCM_SWA_REASON,
    "Phi3VForCausalLM": (
        "ROCm Triton flash attention may run into compilation errors due to "
        "excessive use of shared memory. If this happens, disable Triton FA "
        "by setting `VLLM_USE_TRITON_FLASH_ATTN=0`"
    ),
67
}
68
_ROCM_DEVICE_ID_NAME_MAP: dict[str, str] = {
69
70
71
72
73
74
75
76
    "0x74a0": "AMD_Instinct_MI300A",
    "0x74a1": "AMD_Instinct_MI300X",
    "0x74b5": "AMD_Instinct_MI300X",  # MI300X VF
    "0x74a5": "AMD_Instinct_MI325X",
    "0x74b9": "AMD_Instinct_MI325X",  # MI325X VF
    "0x74a9": "AMD_Instinct_MI300X_HF",
    "0x74bd": "AMD_Instinct_MI300X_HF",
}
77

78
# Prevent use of clashing `{CUDA/HIP}_VISIBLE_DEVICES`
79
80
81
82
83
84
85
86
87
88
89
90
91
92
93
94
95
96
97
98
99
100
101
102
103
if "HIP_VISIBLE_DEVICES" in os.environ:
    val = os.environ["HIP_VISIBLE_DEVICES"]
    if cuda_val := os.environ.get("CUDA_VISIBLE_DEVICES", None):
        assert val == cuda_val
    else:
        os.environ["CUDA_VISIBLE_DEVICES"] = val

# AMDSMI utils
# Note that NVML is not affected by `{CUDA/HIP}_VISIBLE_DEVICES`,
# all the related functions work on real physical device ids.
# the major benefit of using AMDSMI is that it will not initialize CUDA


def with_amdsmi_context(fn):
    @wraps(fn)
    def wrapper(*args, **kwargs):
        amdsmi_init()
        try:
            return fn(*args, **kwargs)
        finally:
            amdsmi_shut_down()

    return wrapper


104
105
106
107
108
109
@cache
def on_gfx1x() -> bool:
    GPU_ARCH = torch.cuda.get_device_properties("cuda").gcnArchName
    return any(arch in GPU_ARCH for arch in ["gfx11", "gfx12"])


110
@cache
111
def on_mi3xx() -> bool:
112
    GPU_ARCH = torch.cuda.get_device_properties("cuda").gcnArchName
113
114
115
116
117
118
119
    return any(arch in GPU_ARCH for arch in ["gfx942", "gfx950"])


@cache
def on_gfx9() -> bool:
    GPU_ARCH = torch.cuda.get_device_properties("cuda").gcnArchName
    return any(arch in GPU_ARCH for arch in ["gfx90a", "gfx942", "gfx950"])
120
121


122
123
124
125
126
127
@cache
def on_gfx950() -> bool:
    GPU_ARCH = torch.cuda.get_device_properties("cuda").gcnArchName
    return any(arch in GPU_ARCH for arch in ["gfx950"])


128
@cache
129
def use_rocm_custom_paged_attention(
130
131
132
133
134
135
136
    qtype: torch.dtype,
    head_size: int,
    block_size: int,
    gqa_ratio: int,
    max_seq_len: int,
    sliding_window: int,
    kv_cache_dtype: str,
137
138
    alibi_slopes: torch.Tensor | None = None,
    sinks: torch.Tensor | None = None,
139
) -> bool:
140
141
    GPU_ARCH = torch.cuda.get_device_properties("cuda").gcnArchName
    ON_GFX9 = any(arch in GPU_ARCH for arch in ["gfx90a", "gfx942", "gfx950"])
142
    ON_GFX11_GFX12 = any(arch in GPU_ARCH for arch in ["gfx11", "gfx12"])
143

144
145
    # custom paged attn always supported on V0. On V1, requires sliding window
    # disabled due to observed numerical discrepancy.
146
    if ON_GFX9:
147
148
149
150
151
152
153
154
155
156
157
        return (
            (not envs.VLLM_USE_V1 or sliding_window == 0 or sliding_window == (-1, -1))
            and (qtype == torch.half or qtype == torch.bfloat16)
            and (head_size == 64 or head_size == 128)
            and (block_size == 16 or block_size == 32)
            and (gqa_ratio >= 1 and gqa_ratio <= 16)
            and max_seq_len <= 128 * 1024
            and (envs.VLLM_ROCM_CUSTOM_PAGED_ATTN)
            and not (envs.VLLM_ROCM_USE_AITER_PAGED_ATTN and envs.VLLM_ROCM_USE_AITER)
            and sinks is None
        )
158
159

    else:
160
161
162
163
164
165
166
167
168
169
170
171
172
173
174
175
176
        return (
            ON_GFX11_GFX12
            and (
                not envs.VLLM_USE_V1
                or sliding_window == 0
                or sliding_window == (-1, -1)
            )
            and (qtype == torch.half or qtype == torch.bfloat16)
            and head_size == 128
            and block_size == 16
            and (gqa_ratio >= 3 and gqa_ratio <= 16)
            and max_seq_len <= 128 * 1024
            and alibi_slopes is None
            and kv_cache_dtype == "auto"
            and envs.VLLM_ROCM_CUSTOM_PAGED_ATTN
            and sinks is None
        )
177
178


179
180
class RocmPlatform(Platform):
    _enum = PlatformEnum.ROCM
181
    device_name: str = "rocm"
182
    device_type: str = "cuda"
183
    dispatch_key: str = "CUDA"
184
    ray_device_key: str = "GPU"
185
    dist_backend: str = "nccl"
186
187
    # rocm shares the same device control env var as CUDA
    device_control_env_var: str = "CUDA_VISIBLE_DEVICES"
188

189
    supported_quantization: list[str] = [
190
191
192
193
194
195
196
197
198
199
200
        "awq",
        "gptq",
        "fp8",
        "compressed-tensors",
        "fbgemm_fp8",
        "gguf",
        "quark",
        "ptpc_fp8",
        "mxfp4",
        "petit_nvfp4",
        "torchao",
201
    ]
202

203
    @classmethod
204
    def get_vit_attn_backend(cls, head_size: int, dtype: torch.dtype) -> "_Backend":
205
        from vllm.attention.backends.registry import _Backend
206
207

        if envs.VLLM_ROCM_USE_AITER and envs.VLLM_ROCM_USE_AITER_MHA and on_gfx9():
208
209
210
            return _Backend.ROCM_AITER_FA
        if on_gfx9():
            return _Backend.FLASH_ATTN
211
212
        return _Backend.TORCH_SDPA

213
    @classmethod
214
215
216
217
218
219
220
221
222
223
224
225
    def get_attn_backend_cls(
        cls,
        selected_backend,
        head_size,
        dtype,
        kv_cache_dtype,
        block_size,
        use_v1,
        use_mla,
        has_sink,
        use_sparse,
    ) -> str:
226
        from vllm.attention.backends.registry import _Backend
227

228
        if use_sparse:
229
            raise NotImplementedError("Sparse Attention is not supported on ROCm.")
230
        if use_mla:
231
232
233
            if not use_v1:
                raise RuntimeError(
                    "MLA attention backends require the V1 engine. "
234
235
                    "Set VLLM_USE_V1=1 to enable them."
                )
236

237
            from vllm.v1.attention.backends.mla.rocm_aiter_mla import (
238
239
                is_aiter_mla_enabled,
            )
240
241

            if selected_backend is None:
242
243
244
245
246
                selected_backend = (
                    _Backend.ROCM_AITER_MLA
                    if is_aiter_mla_enabled() or block_size == 1
                    else _Backend.TRITON_MLA
                )
247
248
249

            if selected_backend == _Backend.TRITON_MLA:
                if block_size != 1:
250
                    logger.info_once("Using Triton MLA backend on V1 engine.")
251
                    return "vllm.v1.attention.backends.mla.triton_mla.TritonMLABackend"
252
253
                raise ValueError(
                    f" The selected backend, {selected_backend.name},"
254
255
                    f"does not support block size {block_size}."
                )
256
            if selected_backend == _Backend.ROCM_AITER_MLA:
257
                if block_size == 1:
258
                    logger.info("Using AITER MLA backend on V1 engine.")
259
260
261
                    return (
                        "vllm.v1.attention.backends.mla.rocm_aiter_mla.AiterMLABackend"  # noqa: E501
                    )
262
263
                raise ValueError(
                    f" The selected backend, {selected_backend.name},"
264
                    f"does not support block size {block_size}."
265
266
                    "(currently only supports block size 1)"
                )
267
268
            raise ValueError(
                f" The selected backend, {selected_backend.name},"
269
270
                f"is not MLA type while requested for MLA backend."
            )
271

272
        if envs.VLLM_USE_V1:
273
274
275
            if selected_backend == _Backend.FLEX_ATTENTION:
                logger.info("Using FlexAttention backend on V1 engine.")
                return "vllm.v1.attention.backends.flex_attention.FlexAttentionBackend"
276
277
278
279
            if (
                envs.VLLM_ROCM_USE_AITER and envs.VLLM_ROCM_USE_AITER_MHA and on_gfx9()
            ) or selected_backend == _Backend.ROCM_AITER_FA:
                logger.info("Using Aiter Flash Attention backend on V1 engine.")
280
281
282
283
                return (
                    "vllm.v1.attention.backends."
                    "rocm_aiter_fa.AiterFlashAttentionBackend"
                )
284
285
286
287
288
289
290
291
292
293
            if (
                envs.VLLM_ROCM_USE_AITER and envs.VLLM_ROCM_USE_AITER_UNIFIED_ATTENTION
            ) or selected_backend == _Backend.ROCM_AITER_UNIFIED_ATTN:
                logger.info("Using Aiter Unified Attention backend on V1 engine.")
                return (
                    "vllm.v1.attention.backends."
                    "rocm_aiter_unified_attn.RocmAiterUnifiedAttentionBackend"
                )
            if (
                envs.VLLM_V1_USE_PREFILL_DECODE_ATTENTION
294
295
                or selected_backend == _Backend.ROCM_ATTN
            ):
296
297
                # rocm specific backend, with aiter and/or
                #   triton prefix-prefill
298
                logger.info("Using Rocm Attention backend on V1 engine.")
299
                return "vllm.v1.attention.backends.rocm_attn.RocmAttentionBackend"
300
301
302
            # default case, using triton unified attention
            logger.info("Using Triton Attention backend on V1 engine.")
            return "vllm.v1.attention.backends.triton_attn.TritonAttentionBackend"
303
304
        raise RuntimeError(
            "V0 attention backends have been removed. Set VLLM_USE_V1=1 "
305
306
            "to select a supported backend."
        )
307

308
309
310
311
312
313
314
    @classmethod
    def set_device(cls, device: torch.device) -> None:
        """
        Set the device for the current platform.
        """
        torch.cuda.set_device(device)

315
    @classmethod
316
    @lru_cache(maxsize=8)
317
    def get_device_capability(cls, device_id: int = 0) -> DeviceCapability | None:
318
319
        major, minor = torch.cuda.get_device_capability(device_id)
        return DeviceCapability(major=major, minor=minor)
320

321
    @classmethod
322
    @with_amdsmi_context
323
    def is_fully_connected(cls, physical_device_ids: list[int]) -> bool:
324
325
326
        """
        Query if the set of gpus are fully connected by xgmi (1 hop)
        """
327
        handles = [amdsmi_get_processor_handles()[i] for i in physical_device_ids]
328
329
330
331
        for i, handle in enumerate(handles):
            for j, peer_handle in enumerate(handles):
                if i < j:
                    try:
332
                        link_type = amdsmi_topo_get_link_type(handle, peer_handle)
333
334
335
336
                        # type is 2 for XGMI
                        if link_type["hops"] != 1 or link_type["type"] != 2:
                            return False
                    except AmdSmiException as error:
337
                        logger.error("AMD 1 hop XGMI detection failed.", exc_info=error)
338
339
340
                        return False
        return True

341
    @classmethod
342
    @with_amdsmi_context
343
    @lru_cache(maxsize=8)
344
    def get_device_name(cls, device_id: int = 0) -> str:
345
        physical_device_id = cls.device_id_to_physical_device_id(device_id)
346
        handle = amdsmi_get_processor_handles()[physical_device_id]
347
348
349
350
351
        asic_info = amdsmi_get_gpu_asic_info(handle)
        device_name: str = asic_info["device_id"]
        if device_name in _ROCM_DEVICE_ID_NAME_MAP:
            return _ROCM_DEVICE_ID_NAME_MAP[device_name]
        return asic_info["market_name"]
352
353
354
355
356

    @classmethod
    def get_device_total_memory(cls, device_id: int = 0) -> int:
        device_props = torch.cuda.get_device_properties(device_id)
        return device_props.total_memory
357
358

    @classmethod
359
    def check_and_update_config(cls, vllm_config: "VllmConfig") -> None:
360
361
        from vllm.config.compilation import CUDAGraphMode

362
        cache_config = vllm_config.cache_config
363
364
365
366
367
        compilation_config = vllm_config.compilation_config
        parallel_config = vllm_config.parallel_config
        is_eager_execution = compilation_config == CUDAGraphMode.NONE

        use_v1 = envs.VLLM_USE_V1
368
369
370
        use_aiter_rms_norm = (
            envs.VLLM_ROCM_USE_AITER and envs.VLLM_ROCM_USE_AITER_RMSNORM
        )
371

372
373
374
        if cache_config and cache_config.block_size is None:
            cache_config.block_size = 16

375
        if parallel_config.worker_cls == "auto":
376
            parallel_config.worker_cls = "vllm.v1.worker.gpu_worker.Worker"
377
        #  Aiter rms norm perform best when CUDA Graph capture is enabled.
378
379
380
381
382
383
        if (
            use_v1
            and use_aiter_rms_norm
            and not is_eager_execution
            and "-rms_norm" not in compilation_config.custom_ops
        ):
384
            compilation_config.custom_ops.append("+rms_norm")
385

386
387
388
    @classmethod
    def verify_model_arch(cls, model_arch: str) -> None:
        if model_arch in _ROCM_UNSUPPORTED_MODELS:
389
390
391
            raise ValueError(
                f"Model architecture '{model_arch}' is not supported by ROCm for now."
            )
392
393
394
395

        if model_arch in _ROCM_PARTIALLY_SUPPORTED_MODELS:
            msg = _ROCM_PARTIALLY_SUPPORTED_MODELS[model_arch]
            logger.warning(
396
397
398
399
                "Model architecture '%s' is partially supported by ROCm: %s",
                model_arch,
                msg,
            )
400

401
402
403
404
405
406
    @classmethod
    def verify_quantization(cls, quant: str) -> None:
        super().verify_quantization(quant)
        if quant == "awq" and not envs.VLLM_USE_TRITON_AWQ:
            logger.warning(
                "Using AWQ quantization with ROCm, but VLLM_USE_TRITON_AWQ"
407
408
                " is not set, enabling VLLM_USE_TRITON_AWQ."
            )
409
        envs.VLLM_USE_TRITON_AWQ = True
410
411
412
413

    @classmethod
    def get_punica_wrapper(cls) -> str:
        return "vllm.lora.punica_wrapper.punica_gpu.PunicaWrapperGPU"
414
415

    @classmethod
416
    def get_current_memory_usage(
417
        cls, device: torch.types.Device | None = None
418
    ) -> float:
419
        torch.cuda.reset_peak_memory_stats(device)
420
        return torch.cuda.mem_get_info(device)[1] - torch.cuda.mem_get_info(device)[0]
421
422
423

    @classmethod
    def get_device_communicator_cls(cls) -> str:
424
425
426
        return (
            "vllm.distributed.device_communicators.cuda_communicator.CudaCommunicator"  # noqa
        )
427

428
429
430
431
432
    @classmethod
    def supports_mx(cls) -> bool:
        gcn_arch = torch.cuda.get_device_properties(0).gcnArchName
        return any(gfx in gcn_arch for gfx in ["gfx95"])

433
434
435
    @classmethod
    def supports_fp8(cls) -> bool:
        gcn_arch = torch.cuda.get_device_properties(0).gcnArchName
436
        return any(gfx in gcn_arch for gfx in ["gfx94", "gfx95", "gfx12"])
437
438
439
440

    @classmethod
    def is_fp8_fnuz(cls) -> bool:
        # only device 0 is checked, this assumes MI300 platforms are homogeneous
441
        return "gfx94" in torch.cuda.get_device_properties(0).gcnArchName
442
443
444
445
446
447
448

    @classmethod
    def fp8_dtype(cls) -> torch.dtype:
        if cls.is_fp8_fnuz():
            return torch.float8_e4m3fnuz
        else:
            return torch.float8_e4m3fn
449

450
451
452
453
    @classmethod
    def use_custom_allreduce(cls) -> bool:
        # We only enable custom allreduce for MI300 series
        gcn_arch = torch.cuda.get_device_properties(0).gcnArchName
454
        supported_archs = ["gfx94", "gfx95"]
455
        return any(gfx in gcn_arch for gfx in supported_archs)
456

457
458
459
460
    @classmethod
    def opaque_attention_op(cls) -> bool:
        return True

461
462
    @classmethod
    def get_cu_count(cls, device_id: int = 0) -> int:
463
        return torch.cuda.get_device_properties(device_id).multi_processor_count
464
465
466

    @classmethod
    def is_navi(cls) -> bool:
467
        return "gfx1" in torch.cuda.get_device_properties(0).gcnArchName
468
469

    @classmethod
470
471
    def get_static_graph_wrapper_cls(cls) -> str:
        return "vllm.compilation.cuda_graph.CUDAGraphWrapper"
472

473
474
475
    @classmethod
    def device_count(cls) -> int:
        return cuda_device_count_stateless()
476
477

    @classmethod
478
479
480
    def is_kv_cache_dtype_supported(
        cls, kv_cache_dtype: str, model_config: "ModelConfig"
    ) -> bool:
481
        return True
482
483

    @classmethod
484
485
    def check_if_supports_dtype(cls, dtype: torch.dtype):
        if dtype == torch.bfloat16:  # noqa: SIM102
486
487
488
489
490
491
492
493
494
495
496
497
498
499
500
            if not cls.has_device_capability(80):
                capability = cls.get_device_capability()
                gpu_name = cls.get_device_name()

                if capability is None:
                    compute_str = "does not have a compute capability"
                else:
                    version_str = capability.as_version_str()
                    compute_str = f"has compute capability {version_str}"

                raise ValueError(
                    "Bfloat16 is only supported on GPUs "
                    "with compute capability of at least 8.0. "
                    f"Your {gpu_name} GPU {compute_str}. "
                    "You can use float16 instead by explicitly setting the "
501
502
                    "`dtype` flag in CLI, for example: --dtype=half."
                )
503
504
505
506

    @classmethod
    def support_hybrid_kv_cache(cls) -> bool:
        return True
507
508
509
510

    @classmethod
    def support_static_graph_mode(cls) -> bool:
        return True