CUDA: fix FA out-of-bounds reads (#7479)
This commit is contained in:
parent
1e374365d1
commit
cd93a28cb1
4 changed files with 8 additions and 8 deletions
|
@ -79,7 +79,7 @@ static __global__ void flash_attn_tile_ext_f32(
|
|||
|
||||
#pragma unroll
|
||||
for (int i0 = 0; i0 < D; i0 += 2*WARP_SIZE) {
|
||||
float2 tmp = Q_f2[j*(nb01/sizeof(float2)) + i0/2 + threadIdx.x];
|
||||
float2 tmp = ic0 + j < ne01 ? Q_f2[j*(nb01/sizeof(float2)) + i0/2 + threadIdx.x] : make_float2(0.0f, 0.0f);
|
||||
Q_f[j][i0 + 0*WARP_SIZE + threadIdx.x] = tmp.x * scale;
|
||||
Q_f[j][i0 + 1*WARP_SIZE + threadIdx.x] = tmp.y * scale;
|
||||
}
|
||||
|
|
Loading…
Add table
Add a link
Reference in a new issue