Skip to content

Commit 43ff5f3

Browse files
committed
mtgpu: disable flash attention on qy1 (MTT S80); disable q3_k and mul_mat_batched_cublas
Signed-off-by: Xiaodong Ye <[email protected]>
1 parent e40b33d commit 43ff5f3

File tree

3 files changed

+25
-1
lines changed

3 files changed

+25
-1
lines changed

ggml/src/ggml-cuda.cu

Lines changed: 14 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -2829,6 +2829,12 @@ GGML_CALL static bool ggml_backend_cuda_supports_op(ggml_backend_t backend, cons
28292829
if (op->op == GGML_OP_MUL_MAT && a->ne[3] != b->ne[3]) {
28302830
return false;
28312831
}
2832+
#ifdef GGML_USE_MUSA
2833+
if (b->type == GGML_TYPE_F16 && b->ne[2]*b->ne[3] > 1 &&
2834+
!ggml_is_transposed(a) && !ggml_is_transposed(b)) {
2835+
return false;
2836+
}
2837+
#endif // GGML_USE_MUSA
28322838
switch (a->type) {
28332839
case GGML_TYPE_F32:
28342840
case GGML_TYPE_F16:
@@ -2852,6 +2858,11 @@ GGML_CALL static bool ggml_backend_cuda_supports_op(ggml_backend_t backend, cons
28522858
case GGML_TYPE_IQ3_XXS:
28532859
case GGML_TYPE_IQ4_NL:
28542860
case GGML_TYPE_IQ4_XS:
2861+
#ifdef GGML_USE_MUSA
2862+
if (a->type == GGML_TYPE_Q3_K) {
2863+
return false;
2864+
}
2865+
#endif // GGML_USE_MUSA
28552866
return true;
28562867
default:
28572868
return false;
@@ -2977,6 +2988,9 @@ GGML_CALL static bool ggml_backend_cuda_supports_op(ggml_backend_t backend, cons
29772988
case GGML_OP_RWKV_WKV:
29782989
return true;
29792990
case GGML_OP_FLASH_ATTN_EXT: {
2991+
#ifndef FLASH_ATTN_AVAILABLE
2992+
return false;
2993+
#endif
29802994
if (op->src[0]->ne[0] == 64 && op->src[1]->type == GGML_TYPE_F16) {
29812995
return true;
29822996
}

ggml/src/ggml-cuda/common.cuh

Lines changed: 6 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -50,6 +50,8 @@
5050
#define CC_RDNA1 (CC_OFFSET_AMD + 1010)
5151
#define CC_RDNA2 (CC_OFFSET_AMD + 1030)
5252
#define CC_RDNA3 (CC_OFFSET_AMD + 1100)
53+
#define CC_QY1 210
54+
#define CC_QY2 220
5355

5456
#define MATRIX_ROW_PADDING 512 // last row of quant. matrices is a multiple of this to avoid out-of-bounds memory accesses
5557

@@ -134,6 +136,10 @@ typedef float2 dfloat2;
134136
#define INT8_MMA_AVAILABLE
135137
#endif // !(defined(GGML_USE_HIPBLAS) && defined(__HIP_PLATFORM_AMD__)) && __CUDA_ARCH__ >= CC_TURING
136138

139+
#if !(defined(GGML_USE_MUSA) && __MUSA_ARCH__ <= CC_QY1)
140+
#define FLASH_ATTN_AVAILABLE
141+
#endif // !(defined(GGML_USE_MUSA) && __MUSA_ARCH__ <= CC_QY1)
142+
137143
static constexpr bool fast_fp16_available(const int cc) {
138144
return cc >= CC_PASCAL && cc != 610;
139145
}

ggml/src/ggml-cuda/fattn-tile-f32.cu

Lines changed: 5 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -44,13 +44,17 @@ static __global__ void flash_attn_tile_ext_f32(
4444
const int ne1,
4545
const int ne2,
4646
const int ne3) {
47+
#ifndef FLASH_ATTN_AVAILABLE
48+
NO_DEVICE_CODE;
49+
return;
50+
#endif // FLASH_ATTN_AVAILABLE
4751
// Skip unused kernel variants for faster compilation:
4852
if (use_logit_softcap && !(D == 128 || D == 256)) {
4953
NO_DEVICE_CODE;
5054
return;
5155
}
5256

53-
//In this kernel Q, K, V are matrices while i, j, k are matrix indices.
57+
// In this kernel Q, K, V are matrices while i, j, k are matrix indices.
5458

5559
const int ic0 = (blockIdx.x / parallel_blocks) * ncols; // Index of the Q/QKV column to work on.
5660
const int ip = blockIdx.x % parallel_blocks; // Index in group of blocks running for the same column in parallel.

0 commit comments

Comments
 (0)