mirror of
https://github.com/ggerganov/llama.cpp.git
synced 2024-12-27 06:39:25 +01:00
cuda : fix FA Q src index (1 -> 0) (#9374)
This commit is contained in:
parent
3f7ccfd649
commit
e079bffb66
@ -152,7 +152,7 @@ static void ggml_cuda_flash_attn_ext_wmma_f16(ggml_backend_cuda_context & ctx, g
|
|||||||
} \
|
} \
|
||||||
|
|
||||||
static void ggml_cuda_flash_attn_ext_vec_f16(ggml_backend_cuda_context & ctx, ggml_tensor * dst) {
|
static void ggml_cuda_flash_attn_ext_vec_f16(ggml_backend_cuda_context & ctx, ggml_tensor * dst) {
|
||||||
ggml_tensor * Q = dst->src[1];
|
ggml_tensor * Q = dst->src[0];
|
||||||
ggml_tensor * K = dst->src[1];
|
ggml_tensor * K = dst->src[1];
|
||||||
ggml_tensor * V = dst->src[2];
|
ggml_tensor * V = dst->src[2];
|
||||||
|
|
||||||
@ -227,7 +227,7 @@ static void ggml_cuda_flash_attn_ext_vec_f16(ggml_backend_cuda_context & ctx, gg
|
|||||||
} \
|
} \
|
||||||
|
|
||||||
static void ggml_cuda_flash_attn_ext_vec_f32(ggml_backend_cuda_context & ctx, ggml_tensor * dst) {
|
static void ggml_cuda_flash_attn_ext_vec_f32(ggml_backend_cuda_context & ctx, ggml_tensor * dst) {
|
||||||
ggml_tensor * Q = dst->src[1];
|
ggml_tensor * Q = dst->src[0];
|
||||||
ggml_tensor * K = dst->src[1];
|
ggml_tensor * K = dst->src[1];
|
||||||
ggml_tensor * V = dst->src[2];
|
ggml_tensor * V = dst->src[2];
|
||||||
|
|
||||||
|
Loading…
Reference in New Issue
Block a user